Skip to content
New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Fix mean subtraction in layer norm kernels #519

Draft
wants to merge 1 commit into
base: main
Choose a base branch
from
Draft
Changes from all commits
Commits
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
8 changes: 5 additions & 3 deletions src/liger_kernel/ops/layer_norm.py
Original file line number Diff line number Diff line change
Expand Up @@ -57,13 +57,14 @@ def _layer_norm_forward_kernel(
B_row = tl.load(B_ptr + col_offsets, mask=mask, other=0)

mean = tl.sum(X_row, axis=0) / n_cols
var = tl.sum((X_row - mean) * (X_row - mean), axis=0) / n_cols
Xmm = tl.where(mask, X_row - mean, 0)
var = tl.sum(Xmm * Xmm, axis=0) / n_cols
rstd = rsqrt(var + eps)

tl.store(Mean_ptr, mean)
tl.store(RSTD_ptr, rstd)

Y_row = (X_row - mean) * rstd * W_row + B_row
Y_row = Xmm * rstd * W_row + B_row

tl.store(Y_ptr + col_offsets, Y_row, mask=mask)

Expand Down Expand Up @@ -118,7 +119,8 @@ def _layer_norm_backward_kernel(
mean = tl.load(Mean_ptr)
rstd = tl.load(RSTD_ptr)

x_hat = (x - mean) * rstd
xmm = tl.where(mask, x - mean, 0)
x_hat = xmm * rstd
wdy = w * dy
c1 = tl.sum(x_hat * wdy, axis=0) / n_cols
c2 = tl.sum(wdy, axis=0) / n_cols
Expand Down