Skip to content
Open
Changes from all commits
Commits
Show all changes
67 commits
Select commit Hold shift + click to select a range
5465d8b
Replace chunked FLA with recurrent gated delta rule for T=1 decode
Gasoonjia Apr 2, 2026
a6ebe8a
Runtime dispatch: recurrent (T=1) vs chunked (T>1) inside triton_op
Gasoonjia Apr 3, 2026
fc5018e
Revert model.py, export.py, main.cpp to main branch
Gasoonjia Apr 3, 2026
c90a8e8
Add tests for recurrent (T=1) and multi-T dispatch
Gasoonjia Apr 3, 2026
ce3e9ca
lint fix - 2
Gasoonjia Apr 3, 2026
8d35c65
lint fix - 2
Gasoonjia Apr 3, 2026
709deb0
Merge branch 'main' into recurrent-fla
Gasoonjia Apr 3, 2026
eff976d
lint fix - 3
Gasoonjia Apr 3, 2026
7dd4280
Optimize recurrent kernel: parallelize over V tiles
Gasoonjia Apr 3, 2026
3a1ee31
Dual-method PTE with GPU-resident state for Qwen3.5 MoE
Apr 5, 2026
63c162e
Use share_mutable_buffers to eliminate select_scatter overhead
Apr 6, 2026
47d6b98
Merge branch 'main' into recurrent-fla
Gasoonjia Apr 6, 2026
375e5c0
lint
Gasoonjia Apr 6, 2026
2b36797
remove reduntdant updates
Gasoonjia Apr 6, 2026
c06d58b
Cross-method AOTI constant sharing for KV cache
Apr 7, 2026
6945b2a
Fix cross-method AOTI constant sharing and add dual-method runner
Gasoonjia Apr 7, 2026
ea51d0d
Remove debug printf and decode_only flag
Gasoonjia Apr 7, 2026
a0a62f1
Lint formatting fixes
Gasoonjia Apr 7, 2026
ca69871
Improve CUDA backend error handling and add dual-method runner fallback
Apr 9, 2026
7c148f7
Add CUDA graph capture/replay for decode method
Apr 10, 2026
ee75c2e
Merge branch 'main' into cuda-graph
Gasoonjia Apr 10, 2026
10e7aad
lint and reformat
Gasoonjia Apr 13, 2026
9042f36
Merge branch 'main' into cuda-graph
Gasoonjia Apr 13, 2026
c19d43e
Add fused GatedDeltaNet decode Triton kernel
Gasoonjia Apr 14, 2026
84d1587
Merge branch 'main' into cuda-graph
Gasoonjia Apr 15, 2026
1c73738
Merge branch 'cuda-graph' into fused-deltanet-decode
Gasoonjia Apr 15, 2026
e00a499
solve claude
Gasoonjia Apr 15, 2026
aa7bb82
Merge branch 'main' into cuda-graph
Gasoonjia Apr 15, 2026
deb1c34
Merge branch 'cuda-graph' into fused-deltanet-decode
Gasoonjia Apr 15, 2026
cef386b
Merge branch 'main' into cuda-graph
Gasoonjia Apr 15, 2026
484ad49
Merge branch 'cuda-graph' into fused-deltanet-decode
Gasoonjia Apr 16, 2026
07be9ee
optimized by KA
Gasoonjia Apr 16, 2026
a342209
lint
Gasoonjia Apr 16, 2026
2d32422
Merge branch 'main' into cuda-graph
Gasoonjia Apr 16, 2026
1270870
Merge branch 'main' into cuda-graph
Gasoonjia Apr 16, 2026
8fc7355
solve stride out of scope
Gasoonjia Apr 17, 2026
2c46ed2
Merge branch 'main' into cuda-graph
Gasoonjia Apr 21, 2026
855eb93
Merge branch 'main' into cuda-graph
Gasoonjia Apr 22, 2026
4237d17
remove unused env var
Gasoonjia Apr 22, 2026
9b4705e
Merge branch 'main' into cuda-graph
Gasoonjia Apr 23, 2026
0492e8d
Add GPU-side Gumbel-max sampling for CUDA graph compatibility
Apr 13, 2026
8c0bbf3
lintrunner
Gasoonjia Apr 13, 2026
5245f64
remove git info
Gasoonjia Apr 23, 2026
63b2ceb
Implement FlashDecoding++ async softmax for split-K SDPA
Gasoonjia Apr 14, 2026
fe64a43
remove git msg
Gasoonjia Apr 15, 2026
c6a4b38
remove tmp files
Gasoonjia Apr 16, 2026
c93f8ae
Revert non-SDPA changes to match main
Apr 16, 2026
39589ae
finetuned using KA
Gasoonjia Apr 16, 2026
1a79d9d
revert KA optimization
Gasoonjia Apr 23, 2026
2ca1b22
Merge branch 'cuda-graph' into fused-deltanet-decode
Gasoonjia Apr 23, 2026
5535d78
Merge branch 'gasoonjia/flashdecoding-pp-async-softmax' into fused-de…
Gasoonjia Apr 23, 2026
ede6d13
try to use KA optimize moe kernel
Gasoonjia Apr 16, 2026
a19ee33
Merge branch 'fused-deltanet-decode' into kamoe
Gasoonjia Apr 23, 2026
880391d
reintro llm headers
Gasoonjia Apr 23, 2026
6f411af
lint
Gasoonjia Apr 24, 2026
eff4294
add top-p and top-k arg
Gasoonjia Apr 24, 2026
61d47aa
move top-p and top-k suport into a individual PR
Gasoonjia Apr 24, 2026
2828ba9
Merge branch 'cuda-graph-sampling' into gasoonjia/flashdecoding-pp-as…
Gasoonjia Apr 24, 2026
f380b22
Merge branch 'gasoonjia/flashdecoding-pp-async-softmax' into fused-de…
Gasoonjia Apr 24, 2026
2dd7810
Merge branch 'fused-deltanet-decode' into kamoe
Gasoonjia Apr 24, 2026
3e185c0
Merge branch 'main' into cuda-graph-sampling
Gasoonjia Apr 27, 2026
e7deb42
Merge branch 'cuda-graph-sampling' into gasoonjia/flashdecoding-pp-as…
Gasoonjia Apr 27, 2026
b263c07
Merge branch 'gasoonjia/flashdecoding-pp-async-softmax' into fused-de…
Gasoonjia Apr 27, 2026
9e1e159
Merge branch 'fused-deltanet-decode' into kamoe
Gasoonjia Apr 27, 2026
4961897
Merge branch 'main' into kamoe
Gasoonjia Apr 27, 2026
dc52e83
Merge branch 'main' into kamoe
Gasoonjia Apr 27, 2026
05cd739
Merge branch 'main' into kamoe
Gasoonjia Apr 27, 2026
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
28 changes: 18 additions & 10 deletions backends/cuda/triton/kernels/fused_moe.py
Original file line number Diff line number Diff line change
Expand Up @@ -51,6 +51,8 @@
triton.Config({"BLOCK_SIZE_N": 8, "BLOCK_SIZE_K": 256}, num_warps=2, num_stages=5),
triton.Config({"BLOCK_SIZE_N": 8, "BLOCK_SIZE_K": 256}, num_warps=2, num_stages=3),
triton.Config({"BLOCK_SIZE_N": 16, "BLOCK_SIZE_K": 256}, num_warps=2, num_stages=5),
triton.Config({"BLOCK_SIZE_N": 16, "BLOCK_SIZE_K": 128}, num_warps=4, num_stages=4),
triton.Config({"BLOCK_SIZE_N": 32, "BLOCK_SIZE_K": 128}, num_warps=4, num_stages=3),
]

# Autotune configs for GEMM2 (_fused_moe_silu_kernel).
Expand All @@ -63,6 +65,8 @@
triton.Config({"BLOCK_SIZE_N": 16, "BLOCK_SIZE_K": 256}, num_warps=4, num_stages=4),
triton.Config({"BLOCK_SIZE_N": 8, "BLOCK_SIZE_K": 256}, num_warps=2, num_stages=3),
triton.Config({"BLOCK_SIZE_N": 8, "BLOCK_SIZE_K": 256}, num_warps=4, num_stages=3),
triton.Config({"BLOCK_SIZE_N": 16, "BLOCK_SIZE_K": 128}, num_warps=2, num_stages=3),
triton.Config({"BLOCK_SIZE_N": 32, "BLOCK_SIZE_K": 128}, num_warps=4, num_stages=4),
]


Expand Down Expand Up @@ -171,9 +175,9 @@ def _fused_moe_kernel(
scale_ptrs, mask=k_mask[:, None] & n_mask[None, :], other=0.0
).to(tl.float32)

# Dequantize and accumulate: vector-matrix multiply
b_dequant = ((b.to(tl.float32) - 8.0) * b_scale).to(compute_type)
acc += tl.sum(a[:, None].to(compute_type) * b_dequant, axis=0)
# Dequantize and accumulate in float32: vector-matrix multiply
b_dequant = (b.to(tl.float32) - 8.0) * b_scale
acc += tl.sum(a[:, None].to(tl.float32) * b_dequant, axis=0)

# Advance K pointers
a_ptrs += BLOCK_SIZE_K * stride_ak
Expand Down Expand Up @@ -259,10 +263,10 @@ def _fused_moe_silu_kernel(
k_remaining = K - k_step * BLOCK_SIZE_K
k_mask = offs_k < k_remaining

# Load gate and up, apply SiLU(gate) * up
# Load gate and up in float32, apply SiLU(gate) * up
gate = tl.load(a_gate_ptrs, mask=k_mask, other=0.0).to(tl.float32)
up = tl.load(a_up_ptrs, mask=k_mask, other=0.0)
a = (gate * tl.sigmoid(gate) * up).to(compute_type)
up = tl.load(a_up_ptrs, mask=k_mask, other=0.0).to(tl.float32)
a = gate * tl.sigmoid(gate) * up

# Load and dequantize INT4 weights
b = tl.load(b_ptrs, mask=k_mask[:, None] & n_mask[None, :], other=0)
Expand Down Expand Up @@ -290,8 +294,8 @@ def _fused_moe_silu_kernel(
scale_ptrs, mask=k_mask[:, None] & n_mask[None, :], other=0.0
).to(tl.float32)

b_dequant = ((b.to(tl.float32) - 8.0) * b_scale).to(compute_type)
acc += tl.sum(a[:, None].to(compute_type) * b_dequant, axis=0)
b_dequant = (b.to(tl.float32) - 8.0) * b_scale
acc += tl.sum(a[:, None] * b_dequant, axis=0)

a_gate_ptrs += BLOCK_SIZE_K * stride_ak
a_up_ptrs += BLOCK_SIZE_K * stride_ak
Expand Down Expand Up @@ -571,6 +575,8 @@ def moe_align_block_size(
triton.Config(
{"BLOCK_SIZE_N": 128, "BLOCK_SIZE_K": 128}, num_warps=4, num_stages=2
),
triton.Config({"BLOCK_SIZE_N": 64, "BLOCK_SIZE_K": 128}, num_warps=8, num_stages=4),
triton.Config({"BLOCK_SIZE_N": 128, "BLOCK_SIZE_K": 64}, num_warps=8, num_stages=4),
]

# Autotune configs for batched GEMM2 (down projection + SiLU).
Expand All @@ -581,6 +587,8 @@ def moe_align_block_size(
triton.Config(
{"BLOCK_SIZE_N": 128, "BLOCK_SIZE_K": 128}, num_warps=4, num_stages=2
),
triton.Config({"BLOCK_SIZE_N": 64, "BLOCK_SIZE_K": 128}, num_warps=8, num_stages=3),
triton.Config({"BLOCK_SIZE_N": 128, "BLOCK_SIZE_K": 64}, num_warps=8, num_stages=4),
]


Expand Down Expand Up @@ -778,9 +786,9 @@ def _fused_moe_silu_batched_kernel(
k_remaining = K - k_step * BLOCK_SIZE_K
k_mask = offs_k < k_remaining

# Load gate and up tiles [BLOCK_M, BLOCK_K], apply SiLU
# Load gate and up in float32 [BLOCK_M, BLOCK_K], apply SiLU
gate = tl.load(a_gate_ptrs, mask=k_mask[None, :], other=0.0).to(tl.float32)
up = tl.load(a_up_ptrs, mask=k_mask[None, :], other=0.0)
up = tl.load(a_up_ptrs, mask=k_mask[None, :], other=0.0).to(tl.float32)
a = (gate * tl.sigmoid(gate) * up).to(compute_type)

# Load and dequantize INT4 weights [BLOCK_K, BLOCK_N]
Expand Down
Loading