@@ -140,17 +140,195 @@ def add_fn(x, y):
140
140
print (f"Vector addition of\n X:\t { x } \n Y:\t { y } \n is equal to\n { out } " )
141
141
142
142
######################################################################
143
- # Composibility and Limitations
143
+ # Composability
144
+ # -------------------------------------------------------------------
145
+ #
146
+ # User-defined triton kernels do not automatically support all PyTorch
147
+ # subsystems, like in the following use cases:
148
+ # - Adding a CPU fallback
149
+ # - Adding a ``FlopCounter`` formula
150
+ # - Composing with Tensor Subclasses
151
+ #
152
+ # To compose with additional PyTorch subsystems, use ``torch.library.triton_op``.
153
+ #
154
+ # triton_op is a structured way of defining a custom operator that is backed by one
155
+ # or more triton kernels: like regular custom operators (``torch.library.custom_op``),
156
+ # you are able to specify the interactions with PyTorch subsystems via ``torch.library``.
157
+ # However, unlike ``torch.library.custom_op``, which creates opaque callables w.r.t.
158
+ # ``torch.compile``, ``torch.compile`` traces into ``triton_op`` to apply optimizations.
159
+ #
160
+ # Here’s a chart of which API to use when integrating triton kernels with PyTorch.
161
+ #
162
+ # .. list-table::
163
+ # :header-rows: 1
164
+ #
165
+ # * -
166
+ # - triton kernel (no explicit torch.library wrapper)
167
+ # - torch.library.triton_op
168
+ # - torch.library.custom_op
169
+ # * - Supports inference
170
+ # - Yes
171
+ # - Yes
172
+ # - Yes
173
+ # * - Supports training
174
+ # - In the majority of cases
175
+ # - Yes
176
+ # - Yes
177
+ # * - Supports torch.compile
178
+ # - Yes
179
+ # - Yes
180
+ # - Yes
181
+ # * - Supports torch.compile(fullgraph=True)
182
+ # - In the majority of cases
183
+ # - In the majority of cases
184
+ # - In all cases
185
+ # * - Does torch.compile trace into the implementation?
186
+ # - Yes
187
+ # - Yes
188
+ # - No
189
+ # * - Supports AOTInductor
190
+ # - Yes
191
+ # - Yes
192
+ # - No
193
+ # * - Supports PyTorch Subsystems like FlopCounterMode, CPU Fallback, Tensor Subclasses
194
+ # - No
195
+ # - Yes
196
+ # - Yes
197
+
198
+ ######################################################################
199
+ # Wrapping triton kernels with triton_op
200
+ # ^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^^
201
+ #
202
+ # Use ``torch.library.triton_op`` to wrap a function that may invoke one or more triton kernels.
203
+ # Use ``torch.library.wrap_triton`` to wrap the calls to the triton kernel.
204
+
205
+ from torch .library import triton_op , wrap_triton
206
+
207
+ @triton_op ("mylib::mysin" , mutates_args = {})
208
+ def mysin (x ):
209
+ out = torch .empty_like (x )
210
+ n_elements = x .numel ()
211
+ wrap_triton (sin_kernel )[(n_elements ,)](x , out , n_elements , BLOCK_SIZE = 4 )
212
+ return out
213
+
214
+ ######################################################################
215
+ # You can invoke the ``triton_op`` in one of the following two ways.
216
+
217
+ x = torch .randn (3 , device = "cuda" )
218
+ y = mysin (x )
219
+ z = torch .ops .mylib .mysin .default (x )
220
+
221
+ assert torch .allclose (y , x .sin ())
222
+ assert torch .allclose (z , x .sin ())
223
+
224
+ ######################################################################
225
+ # The resulting ``triton_op`` works with ``torch.compile`` and ``AOTInductor``.
226
+
227
+ y = torch .compile (mysin )(x )
228
+ assert torch .allclose (y , x .sin ())
229
+
230
+ ######################################################################
231
+ # Adding training support
232
+ # ^^^^^^^^^^^^^^^^^^^^^^^
233
+ #
234
+ # Use ``register_autograd`` to add an autograd formula for the ``triton_op``.
235
+ # Prefer this to using ``torch.autograd.Function`` (which has various composability footguns
236
+ # with ``torch.compile``).
237
+
238
+ def backward (ctx , grad_output ):
239
+ x , = ctx .saved_tensors
240
+ return grad_input * x .cos ()
241
+
242
+ def setup_context (ctx , inputs , output ):
243
+ x , = inputs
244
+ ctx .save_for_backward (x )
245
+
246
+ mysin .register_autograd (backward , setup_context = setup_context )
247
+
248
+ ######################################################################
249
+ # Note that the backward must be a composition of PyTorch-understood operators.
250
+ # If you want the backward to call triton kernels, then those must be wrapped in ``triton_op`` as well:
251
+
252
+ @triton .jit
253
+ def cos_kernel (
254
+ in_ptr0 ,
255
+ out_ptr ,
256
+ n_elements ,
257
+ BLOCK_SIZE : "tl.constexpr" ,
258
+ ):
259
+ pid = tl .program_id (axis = 0 )
260
+ block_start = pid * BLOCK_SIZE
261
+ offsets = block_start + tl .arange (0 , BLOCK_SIZE )
262
+ mask = offsets < n_elements
263
+ x = tl .load (in_ptr0 + offsets , mask = mask )
264
+ output = tl .cos (x )
265
+ tl .store (out_ptr + offsets , output , mask = mask )
266
+
267
+ @triton_op ("mylib::mycos" , mutates_args = {})
268
+ def mycos (x ):
269
+ out = torch .empty_like (x )
270
+ n_elements = x .numel ()
271
+ wrap_triton (cos_kernel )[(n_elements ,)](x , out , n_elements , BLOCK_SIZE = 4 )
272
+ return out
273
+
274
+ def backward (ctx , grad_output ):
275
+ x , = ctx .saved_tensors
276
+ return grad_input * mycos (x )
277
+
278
+ def setup_context (ctx , inputs , output ):
279
+ x , = inputs
280
+ ctx .save_for_backward (x )
281
+
282
+ mysin .register_autograd (backward , setup_context = setup_context )
283
+
284
+ ######################################################################
285
+ # Adding a CPU Fallback
286
+ # ^^^^^^^^^^^^^^^^^^^^^
287
+ # triton kernels don’t run on CPU. Use ``register_kernel`` to add a CPU (or any other device) fallback for the ``triton_op``:
288
+
289
+ @mysin .register_kernel ("cpu" )
290
+ def _ (x ):
291
+ return torch .sin (x )
292
+
293
+ x = torch .randn (3 )
294
+ y = mysin (x )
295
+ assert torch .allclose (y , x .sin ())
296
+
297
+ The fallback must be composed of PyTorch operators .
298
+
299
+ ######################################################################
300
+ # Adding a FlopCounter Formula
301
+ # ^^^^^^^^^^^^^^^^^^^^^^^^^^^^
302
+ #
303
+ # To specify how many flops the triton kernel reports under PyTorch's flop counter,
304
+ # use ``register_flop_formula``.
305
+
306
+ from torch .utils .flop_counter import FlopCounterMode , register_flop_formula
307
+
308
+ @register_flop_formula (torch .ops .mylib .mysin )
309
+ def _ (x_shape ):
310
+ numel = 1
311
+ for s in x_shape :
312
+ numel *= s
313
+ return numel
314
+
315
+ x = torch .randn (3 , device = "cuda" )
316
+ with FlopCounterMode () as flop_counter :
317
+ y = mysin (x )
318
+
319
+ ######################################################################
320
+ # Limitations
144
321
# --------------------------------------------------------------------
145
322
#
146
323
# As of PyTorch 2.3, the support for user-defined Triton kernels in ``torch.compile``
147
324
# includes dynamic shapes, ``torch.autograd.Function``, JIT inductor, and AOT inductor.
148
325
# You can use these features together to build complex, high-performance models.
149
326
#
327
+ # PyTorch 2.6 added ``torch.library.triton_op``, which adds support for
328
+ # user-defined Triton kernels in tensor subclasses and other advanced features.
329
+ #
150
330
# However, there are certain limitations to be aware of:
151
331
#
152
- # * **Tensor Subclasses:** Currently, there is no support for
153
- # tensor subclasses and other advanced features.
154
332
# * **Triton Features:** While ``triton.heuristics`` can be used either standalone or
155
333
# before ``triton.autotune``, it cannot be used after ``triton.autotune``. This
156
334
# implies that if ``triton.heuristics`` and ``triton.autotune`` are to be used
0 commit comments