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

[Pallas] Improve FlashAttention segment_ids test case #7034

Merged
merged 6 commits into from
May 8, 2024
Merged
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
33 changes: 16 additions & 17 deletions test/test_pallas.py
Original file line number Diff line number Diff line change
Expand Up @@ -643,24 +643,23 @@ def test_flash_attention_wrapper_segment_ids_1(self):
q = torch.randn(3, 2, 128, 4)
k = torch.randn(3, 2, 128, 4)
v = torch.randn(3, 2, 128, 4)
q_segment_ids = torch.zeros(3, 128)
kv_segment_ids = torch.zeros(3, 128)
zeros = torch.zeros(3, 32)
segment_ids = torch.cat([zeros, zeros + 1, zeros + 2, zeros + 3], dim=1)
o = flash_attention(
q.to("xla"), k.to("xla"), v.to("xla"), False, q_segment_ids.to("xla"),
kv_segment_ids.to("xla"))
q.to("xla"), k.to("xla"), v.to("xla"), False, segment_ids.to("xla"),
segment_ids.to("xla"))

jax_q = jnp.array(q.numpy(), dtype=jnp.float32)
jax_k = jnp.array(k.numpy(), dtype=jnp.float32)
jax_v = jnp.array(v.numpy(), dtype=jnp.float32)
jax_q_segment_ids = jnp.array(q_segment_ids.numpy(), dtype=jnp.float32)
jax_kv_segment_ids = jnp.array(kv_segment_ids.numpy(), dtype=jnp.float32)
jax_segment_ids = jnp.array(segment_ids.numpy(), dtype=jnp.float32)
expected_o = torch.from_numpy(
np.array(
jax_flash_attention(
jax_q,
jax_k,
jax_v,
segment_ids=SegmentIds(jax_q_segment_ids, jax_kv_segment_ids),
segment_ids=SegmentIds(jax_segment_ids, jax_segment_ids),
)))

self.assertTrue(torch.allclose(o.cpu(), expected_o.cpu(), atol=1e-05))
Expand All @@ -674,16 +673,16 @@ def test_flash_attention_wrapper_segment_ids_2(self):
q = torch.randn(3, 2, 128, 4).to("xla")
k = torch.randn(3, 2, 128, 4).to("xla")
v = torch.randn(3, 2, 128, 4).to("xla")
q_segment_ids = torch.zeros(3, 128).to("xla")
kv_segment_ids = torch.zeros(3, 128).to("xla")
o = flash_attention(q, k, v, False, q_segment_ids, kv_segment_ids)
zeros = torch.zeros(3, 32).to("xla")
segment_ids = torch.cat([zeros, zeros + 1, zeros + 2, zeros + 3], dim=1)
o = flash_attention(q, k, v, False, segment_ids, segment_ids)

expected_o = self._attention(
q,
k,
v,
attn_mask=self._make_attention_mask_from_segment_ids(
q_segment_ids, kv_segment_ids))
segment_ids, segment_ids))
self.assertTrue(torch.allclose(o.cpu(), expected_o.cpu(), atol=1e-05))
jax.config.update('jax_default_matmul_precision', jax.lax.Precision.DEFAULT)

Expand All @@ -697,13 +696,13 @@ def test_flash_attention_backward_segment_ids(self):
q = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
k = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
v = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
q_segment_ids = torch.zeros(4, 128).to("xla")
kv_segment_ids = torch.zeros(4, 128).to("xla")
zeros = torch.zeros(4, 32).to("xla")
segment_ids = torch.cat([zeros, zeros + 1, zeros + 2, zeros + 3], dim=1)
q.retain_grad()
k.retain_grad()
v.retain_grad()

o = flash_attention(q, k, v, False, q_segment_ids, kv_segment_ids)
o = flash_attention(q, k, v, False, segment_ids, segment_ids)
loss = o.sum()
loss.backward()
xm.mark_step()
Expand All @@ -716,8 +715,8 @@ def test_flash_attention_backward_segment_ids(self):
q = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
k = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
v = torch.randn(4, 2, 128, 8, requires_grad=True).to("xla")
q_segment_ids = torch.zeros(4, 128).to("xla")
kv_segment_ids = torch.zeros(4, 128).to("xla")
zeros = torch.zeros(4, 32).to("xla")
segment_ids = torch.cat([zeros, zeros + 1, zeros + 2, zeros + 3], dim=1)
q.retain_grad()
k.retain_grad()
v.retain_grad()
Expand All @@ -727,7 +726,7 @@ def test_flash_attention_backward_segment_ids(self):
k,
v,
attn_mask=self._make_attention_mask_from_segment_ids(
q_segment_ids, kv_segment_ids))
segment_ids, segment_ids))
loss = o.sum()
loss.backward()
xm.mark_step()
Expand Down
Loading