Skip to content

Commit aafe47a

Browse files
committed
[example] fused_linear_jsd and tests
1 parent 1e94f92 commit aafe47a

File tree

5 files changed

+251
-1
lines changed

5 files changed

+251
-1
lines changed

benchmarks/run.py

Lines changed: 6 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -151,6 +151,11 @@ class RunResult:
151151
("examples.matmul_split_k", "matmul_split_k_tritonbench"),
152152
],
153153
),
154+
"fused_linear_jsd": (
155+
"tritonbench.operators.fused_linear_jsd.operator",
156+
"examples.fused_linear_jsd",
157+
"fused_linear_jsd_fwd_tritonbench",
158+
),
154159
}
155160

156161

@@ -519,6 +524,7 @@ def helion_method(
519524
attr.settings.force_autotune = True
520525
attr.settings.static_shape = True # pyright: ignore[reportAttributeAccessIssue]
521526

527+
kfunc._self = self # pyright: ignore[reportFunctionMemberAccess]
522528
if isinstance(kfunc, Kernel):
523529
# Helion kernel - we call it in a lambda to delay execution until measurement
524530
measured_func_callable = lambda: kfunc(*args, **kwargs) # noqa: E731

examples/fused_linear_jsd.py

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

helion/autotuner/base_search.py

Lines changed: 2 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -53,6 +53,7 @@
5353
"misaligned address", # CUDA Error
5454
"PassManager::run failed", # Triton Error
5555
"illegal memory access", # CUDA Error
56+
"exceeds triton maximum tensor numel", # Triton Error
5657
],
5758
)
5859
)
@@ -149,7 +150,7 @@ def benchmark_function(self, config: Config, fn: CompiledConfig) -> float:
149150
self.log.warning(f"PTXASError compiling config: {config}")
150151
except Exception as e:
151152
msg = str(e)
152-
if not _expected_errors_regexp.search(msg):
153+
if not _expected_errors_regexp.search(msg + str(e.__cause__)):
153154
raise exc.TritonError(f"{type(e).__qualname__}: {e}", config) from e
154155
# Surface Triton IR pass failures more prominently for easier bug reports.
155156
if "PassManager::run failed" in msg:

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_jagged_dense_add)
845917
from __future__ import annotations
846918

test/test_examples.py

Lines changed: 35 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1042,6 +1042,41 @@ def test_swiglu(self):
10421042
)
10431043
)
10441044

1045+
def test_fused_linear_jsd(self):
1046+
beta = 0.5
1047+
ignore_index = 1
1048+
temperature = 1.0
1049+
m, n, k = 64, 128, 256
1050+
1051+
student_input = torch.randn([m, n], device=DEVICE, dtype=torch.float32)
1052+
teacher_input = torch.randn([m, n], device=DEVICE, dtype=torch.float32)
1053+
student_weight = torch.randn([k, n], device=DEVICE, dtype=torch.float32)
1054+
teacher_weight = torch.randn([k, n], device=DEVICE, dtype=torch.float32)
1055+
1056+
args = (
1057+
beta,
1058+
ignore_index,
1059+
temperature,
1060+
student_weight,
1061+
teacher_weight,
1062+
student_input,
1063+
teacher_input,
1064+
)
1065+
1066+
# Import and use the reference implementation
1067+
mod = import_path(EXAMPLES_DIR / "fused_linear_jsd.py")
1068+
expected = mod.fused_linear_jsd_pytorch(*args)
1069+
1070+
self.assertExpectedJournal(
1071+
check_example(
1072+
"fused_linear_jsd",
1073+
args,
1074+
expected,
1075+
fn_name="fused_linear_jsd_fwd",
1076+
block_sizes=[32],
1077+
)
1078+
)
1079+
10451080

10461081
if __name__ == "__main__":
10471082
unittest.main()

0 commit comments

Comments
 (0)