Skip to content

Commit 6bd9219

Browse files
committed
[example] fused_linear_jsd and tests
1 parent 8442721 commit 6bd9219

File tree

4 files changed

+251
-0
lines changed

4 files changed

+251
-0
lines changed

benchmarks/run.py

Lines changed: 5 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -180,6 +180,11 @@ class RunResult:
180180
"examples.jagged_softmax",
181181
"jagged_softmax_tritonbench",
182182
),
183+
"fused_linear_jsd": (
184+
"tritonbench.operators.fused_linear_jsd.operator",
185+
"examples.fused_linear_jsd",
186+
"fused_linear_jsd_fwd_tritonbench",
187+
),
183188
# Multiple kernel variants:
184189
"gemm": (
185190
"tritonbench.operators.gemm.operator",

examples/fused_linear_jsd.py

Lines changed: 139 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -0,0 +1,139 @@
1+
"""
2+
Fused Linear JSD Example
3+
===========================
4+
5+
This example demonstrates how to implement a JSD kernel using Helion and
6+
fuse it with a linear layer.
7+
"""
8+
9+
# %%
10+
# Imports
11+
# -------
12+
from __future__ import annotations
13+
14+
from typing import Callable
15+
16+
import torch
17+
18+
import helion
19+
from helion._testing import run_example
20+
import helion.language as hl
21+
22+
23+
# %%
24+
# Helion Kernel
25+
# -------------------
26+
@helion.kernel()
27+
def fused_linear_jsd_fwd(
28+
beta: float,
29+
ignore_index: int,
30+
temperature: float,
31+
student_weight: torch.Tensor,
32+
teacher_weight: torch.Tensor,
33+
student_input: torch.Tensor,
34+
teacher_input: torch.Tensor,
35+
) -> torch.Tensor:
36+
student_logits = student_input @ student_weight.T
37+
teacher_logits = teacher_input @ teacher_weight.T
38+
loss = student_logits.new_empty(student_input.shape[0], dtype=torch.float)
39+
for batch in hl.tile(student_logits.shape[0]):
40+
student_prob = torch.log_softmax(student_logits[batch, :] / temperature, dim=-1)
41+
teacher_prob = torch.log_softmax(teacher_logits[batch, :] / temperature, dim=-1)
42+
student_prob = student_prob.to(torch.float).view(-1, student_prob.size(-1))
43+
teacher_prob = teacher_prob.to(torch.float).view(-1, teacher_prob.size(-1))
44+
m = torch.exp(student_prob) + beta * (
45+
torch.exp(teacher_prob) - torch.exp(student_prob)
46+
)
47+
teacher_div = torch.nn.functional.kl_div(
48+
torch.log(m), teacher_prob, reduction="none", log_target=True
49+
).sum(dim=-1)
50+
student_div = torch.nn.functional.kl_div(
51+
torch.log(m), student_prob, reduction="none", log_target=True
52+
).sum(dim=-1)
53+
batch_loss = student_div + beta * (teacher_div - student_div)
54+
loss[batch] = batch_loss
55+
return (loss / student_logits.shape[0]).sum()
56+
57+
58+
# %%
59+
# Benchmark Entry Point Function
60+
# -------------------
61+
def fused_linear_jsd_fwd_tritonbench(
62+
tb_op: object,
63+
student_input: torch.Tensor,
64+
teacher_input: torch.Tensor,
65+
label: torch.Tensor | None = None,
66+
) -> Callable[[], torch.Tensor]:
67+
assert label is None
68+
baseline_op = tb_op.baseline_op # pyright: ignore[reportAttributeAccessIssue]
69+
beta = baseline_op.jsd.beta
70+
ignore_index = baseline_op.jsd.ignore_index
71+
temperature = baseline_op.temperature
72+
student_weight = baseline_op.student_lin.weight
73+
teacher_weight = baseline_op.teacher_lin.weight
74+
return lambda: fused_linear_jsd_fwd(
75+
beta,
76+
ignore_index,
77+
temperature,
78+
student_weight,
79+
teacher_weight,
80+
student_input,
81+
teacher_input,
82+
)
83+
84+
85+
# %%
86+
# Reference Implementation
87+
# --------------------
88+
def fused_linear_jsd_pytorch(
89+
beta: float,
90+
ignore_index: int,
91+
temperature: float,
92+
student_weight: torch.Tensor,
93+
teacher_weight: torch.Tensor,
94+
student_input: torch.Tensor,
95+
teacher_input: torch.Tensor,
96+
) -> torch.Tensor:
97+
student_logits = student_input @ student_weight.T
98+
teacher_logits = teacher_input @ teacher_weight.T
99+
student_prob = torch.log_softmax(student_logits / temperature, dim=-1)
100+
teacher_prob = torch.log_softmax(teacher_logits / temperature, dim=-1)
101+
student_prob = student_prob.to(torch.float).view(-1, student_prob.size(-1))
102+
teacher_prob = teacher_prob.to(torch.float).view(-1, teacher_prob.size(-1))
103+
m = torch.exp(student_prob) + beta * (
104+
torch.exp(teacher_prob) - torch.exp(student_prob)
105+
)
106+
teacher_div = torch.nn.functional.kl_div(
107+
torch.log(m), teacher_prob, reduction="none", log_target=True
108+
).sum(dim=-1)
109+
student_div = torch.nn.functional.kl_div(
110+
torch.log(m), student_prob, reduction="none", log_target=True
111+
).sum(dim=-1)
112+
loss = student_div + beta * (teacher_div - student_div)
113+
return (loss / student_logits.shape[0]).sum()
114+
115+
116+
# %%
117+
# Verification Function
118+
# -------------------
119+
def check(m: int, n: int, k: int) -> None:
120+
student_input = torch.rand([m, n], device="cuda", dtype=torch.float)
121+
teacher_input = torch.rand([m, n], device="cuda", dtype=torch.float)
122+
student_weight = torch.rand([k, n], device="cuda", dtype=torch.float)
123+
teacher_weight = torch.rand([k, n], device="cuda", dtype=torch.float)
124+
run_example(
125+
fused_linear_jsd_fwd,
126+
fused_linear_jsd_pytorch,
127+
(0.5, 1, 1.0, student_weight, teacher_weight, student_input, teacher_input),
128+
)
129+
130+
131+
# %%
132+
# Main Function
133+
# -----------
134+
def main() -> None:
135+
check(1024, 4096, 128256)
136+
137+
138+
if __name__ == "__main__":
139+
main()

test/test_examples.expected

Lines changed: 72 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -841,6 +841,78 @@ def fp8_gemm(x: torch.Tensor, y: torch.Tensor, *, _launcher=_default_launcher):
841841
_launcher(_helion_fp8_gemm, (triton.cdiv(256, _BLOCK_SIZE_0) * triton.cdiv(256, _BLOCK_SIZE_1),), x, y, out, _BLOCK_SIZE_0, _BLOCK_SIZE_1, _BLOCK_SIZE_2, num_warps=4, num_stages=3)
842842
return out
843843

844+
--- assertExpectedJournal(TestExamples.test_fused_linear_jsd)
845+
from __future__ import annotations
846+
847+
import torch
848+
import triton
849+
import triton.language as tl
850+
from torch._inductor.runtime.triton_helpers import math as tl_math
851+
from torch._inductor.runtime.triton_compat import libdevice
852+
from helion.runtime import default_launcher as _default_launcher
853+
854+
@triton.jit
855+
def _helion_fused_linear_jsd_fwd(student_logits, teacher_logits, loss, student_input_size_0, student_weight_size_0, loss_stride_0, student_logits_stride_0, student_logits_stride_1, teacher_logits_stride_0, teacher_logits_stride_1, temperature, beta, _BLOCK_SIZE_0: tl.constexpr, _RDIM_SIZE_1: tl.constexpr):
856+
pid_0 = tl.program_id(0)
857+
offset_0 = pid_0 * _BLOCK_SIZE_0
858+
indices_0 = (offset_0 + tl.arange(0, _BLOCK_SIZE_0)).to(tl.int32)
859+
mask_0 = indices_0 < student_input_size_0
860+
indices_1 = tl.arange(0, _RDIM_SIZE_1).to(tl.int32)
861+
mask_1 = indices_1 < student_weight_size_0
862+
load = tl.load(student_logits + (indices_0[:, None] * student_logits_stride_0 + indices_1[None, :] * student_logits_stride_1), mask_0[:, None] & mask_1[None, :], other=0)
863+
v_0 = load / temperature
864+
_mask_to = tl.where(mask_0[:, None] & mask_1[None, :], v_0, tl.full([], float('-inf'), tl.float32))
865+
amax = tl.cast(tl.reshape(tl.max(_mask_to, 1), [_BLOCK_SIZE_0, 1]), tl.float32)
866+
v_1 = v_0 - amax
867+
v_2 = libdevice.exp(v_1)
868+
_mask_to_1 = tl.where(mask_0[:, None] & mask_1[None, :], v_2, tl.full([], 0, tl.float32))
869+
sum_1 = tl.cast(tl.reshape(tl.sum(_mask_to_1, 1), [_BLOCK_SIZE_0, 1]), tl.float32)
870+
v_3 = tl_math.log(sum_1)
871+
v_4 = v_1 - v_3
872+
load_1 = tl.load(teacher_logits + (indices_0[:, None] * teacher_logits_stride_0 + indices_1[None, :] * teacher_logits_stride_1), mask_0[:, None] & mask_1[None, :], other=0)
873+
v_5 = load_1 / temperature
874+
_mask_to_2 = tl.where(mask_0[:, None] & mask_1[None, :], v_5, tl.full([], float('-inf'), tl.float32))
875+
amax_1 = tl.cast(tl.reshape(tl.max(_mask_to_2, 1), [_BLOCK_SIZE_0, 1]), tl.float32)
876+
v_6 = v_5 - amax_1
877+
v_7 = libdevice.exp(v_6)
878+
_mask_to_3 = tl.where(mask_0[:, None] & mask_1[None, :], v_7, tl.full([], 0, tl.float32))
879+
sum_2 = tl.cast(tl.reshape(tl.sum(_mask_to_3, 1), [_BLOCK_SIZE_0, 1]), tl.float32)
880+
v_8 = tl_math.log(sum_2)
881+
v_9 = v_6 - v_8
882+
student_prob_1 = tl.reshape(v_4, [_BLOCK_SIZE_0, _RDIM_SIZE_1])
883+
teacher_prob_1 = tl.reshape(v_9, [_BLOCK_SIZE_0, _RDIM_SIZE_1])
884+
v_10 = libdevice.exp(student_prob_1)
885+
v_11 = libdevice.exp(teacher_prob_1)
886+
v_12 = libdevice.exp(student_prob_1)
887+
v_13 = v_11 - v_12
888+
v_14 = v_13 * beta
889+
v_15 = v_10 + v_14
890+
v_16 = tl_math.log(v_15)
891+
v_17 = teacher_prob_1 - v_16
892+
v_18 = libdevice.exp(teacher_prob_1)
893+
v_19 = v_18 * v_17
894+
_mask_to_4 = tl.where(mask_0[:, None] & mask_1[None, :], v_19, tl.full([], 0, tl.float32))
895+
teacher_div = tl.cast(tl.sum(_mask_to_4, 1), tl.float32)
896+
v_20 = tl_math.log(v_15)
897+
v_21 = student_prob_1 - v_20
898+
v_22 = libdevice.exp(student_prob_1)
899+
v_23 = v_22 * v_21
900+
_mask_to_5 = tl.where(mask_0[:, None] & mask_1[None, :], v_23, tl.full([], 0, tl.float32))
901+
student_div = tl.cast(tl.sum(_mask_to_5, 1), tl.float32)
902+
v_24 = teacher_div - student_div
903+
v_25 = v_24 * beta
904+
v_26 = student_div + v_25
905+
tl.store(loss + indices_0 * loss_stride_0, v_26, mask_0)
906+
907+
def fused_linear_jsd_fwd(beta: float, ignore_index: int, temperature: float, student_weight: torch.Tensor, teacher_weight: torch.Tensor, student_input: torch.Tensor, teacher_input: torch.Tensor, *, _launcher=_default_launcher):
908+
student_logits = student_input @ student_weight.T
909+
teacher_logits = teacher_input @ teacher_weight.T
910+
loss = student_logits.new_empty(student_input.shape[0], dtype=torch.float)
911+
_BLOCK_SIZE_0 = 32
912+
_RDIM_SIZE_1 = triton.next_power_of_2(student_weight.size(0))
913+
_launcher(_helion_fused_linear_jsd_fwd, (triton.cdiv(student_input.size(0), _BLOCK_SIZE_0),), student_logits, teacher_logits, loss, student_input.size(0), student_weight.size(0), loss.stride(0), student_logits.stride(0), student_logits.stride(1), teacher_logits.stride(0), teacher_logits.stride(1), temperature, beta, _BLOCK_SIZE_0, _RDIM_SIZE_1, num_warps=4, num_stages=3)
914+
return (loss / student_logits.shape[0]).sum()
915+
844916
--- assertExpectedJournal(TestExamples.test_geglu)
845917
from __future__ import annotations
846918

test/test_examples.py

Lines changed: 35 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1147,6 +1147,41 @@ def test_int4_gemm(self):
11471147
)
11481148
)
11491149

1150+
def test_fused_linear_jsd(self):
1151+
beta = 0.5
1152+
ignore_index = 1
1153+
temperature = 1.0
1154+
m, n, k = 64, 128, 256
1155+
1156+
student_input = torch.randn([m, n], device=DEVICE, dtype=torch.float32)
1157+
teacher_input = torch.randn([m, n], device=DEVICE, dtype=torch.float32)
1158+
student_weight = torch.randn([k, n], device=DEVICE, dtype=torch.float32)
1159+
teacher_weight = torch.randn([k, n], device=DEVICE, dtype=torch.float32)
1160+
1161+
args = (
1162+
beta,
1163+
ignore_index,
1164+
temperature,
1165+
student_weight,
1166+
teacher_weight,
1167+
student_input,
1168+
teacher_input,
1169+
)
1170+
1171+
# Import and use the reference implementation
1172+
mod = import_path(EXAMPLES_DIR / "fused_linear_jsd.py")
1173+
expected = mod.fused_linear_jsd_pytorch(*args)
1174+
1175+
self.assertExpectedJournal(
1176+
check_example(
1177+
"fused_linear_jsd",
1178+
args,
1179+
expected,
1180+
fn_name="fused_linear_jsd_fwd",
1181+
block_sizes=[32],
1182+
)
1183+
)
1184+
11501185

11511186
if __name__ == "__main__":
11521187
unittest.main()

0 commit comments

Comments
 (0)