-
Notifications
You must be signed in to change notification settings - Fork 124
[Train] Support custom attention with context parallel #910
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
Open
lxd-cumt
wants to merge
4
commits into
flagos-ai:main
Choose a base branch
from
lxd-cumt:ring_flex_attn
base: main
Could not load branches
Branch not found: {{ refName }}
Loading
Could not load tags
Nothing to show
Loading
Are you sure you want to change the base?
Some commits from the old base branch may be removed from the timeline,
and old review comments may become outdated.
Open
Changes from all commits
Commits
Show all changes
4 commits
Select commit
Hold shift + click to select a range
File filter
Filter by extension
Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
There are no files selected for viewing
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
213 changes: 213 additions & 0 deletions
213
flagscale/backends/Megatron-LM/megatron/core/transformer/flex_attention.py.patch
This file contains hidden or bidirectional Unicode text that may be interpreted or compiled differently than what appears below. To review, open the file in an editor that reveals hidden Unicode characters.
Learn more about bidirectional Unicode characters
| Original file line number | Diff line number | Diff line change |
|---|---|---|
| @@ -0,0 +1,213 @@ | ||
| diff --git a/megatron/core/transformer/flex_attention.py b/megatron/core/transformer/flex_attention.py | ||
| new file mode 100644 | ||
| index 000000000..268df27d5 | ||
| --- /dev/null | ||
| +++ b/megatron/core/transformer/flex_attention.py | ||
| @@ -0,0 +1,207 @@ | ||
| +# Copyright (c) 2025, BAAI. All rights reserved. | ||
| + | ||
| +from abc import ABC, abstractmethod | ||
| +from dataclasses import dataclass | ||
| +from typing import NoReturn, Optional, Tuple, Union | ||
| + | ||
| +import torch | ||
| +from torch import Tensor | ||
| + | ||
| +from megatron.core import tensor_parallel | ||
| +from megatron.core.inference.contexts import BaseInferenceContext | ||
| +from megatron.core.models.common.embeddings.rope_utils import ( | ||
| + apply_rotary_pos_emb, | ||
| + apply_rotary_pos_emb_with_cos_sin, | ||
| +) | ||
| +from megatron.core.packed_seq_params import PackedSeqParams | ||
| +from megatron.core.parallel_state import ( | ||
| + get_context_parallel_group, | ||
| +) | ||
| +from megatron.core.utils import ( | ||
| + deprecate_inference_params, | ||
| + divide, | ||
| + get_pg_size, | ||
| + is_fa_min_version, | ||
| + is_te_min_version, | ||
| + nvtx_range_pop, | ||
| + nvtx_range_push, | ||
| +) | ||
| +from .enums import AttnMaskType | ||
| +from .transformer_config import TransformerConfig | ||
| +from megatron.core.process_groups_config import ProcessGroupCollection | ||
| +from megatron.core.transformer.attention import SelfAttention | ||
| +from megatron.core.transformer.attention import SelfAttentionSubmodules | ||
| + | ||
| +from megatron.core.transformer.ring_attention import ring_attn | ||
| + | ||
| + | ||
| +class FlexAttention(SelfAttention): | ||
| + def __init__( | ||
| + self, | ||
| + config: TransformerConfig, | ||
| + submodules: SelfAttentionSubmodules, | ||
| + layer_number: int, | ||
| + attn_mask_type=AttnMaskType.padding, | ||
| + cp_comm_type: str = None, | ||
| + pg_collection: ProcessGroupCollection = None, | ||
| + ): | ||
| + super().__init__( | ||
| + config=config, | ||
| + submodules=submodules, | ||
| + layer_number=layer_number, | ||
| + attn_mask_type=attn_mask_type, | ||
| + cp_comm_type=cp_comm_type, | ||
| + pg_collection=pg_collection, | ||
| + ) | ||
| + | ||
| + self.config = config | ||
| + | ||
| + def forward( | ||
| + self, | ||
| + hidden_states: Tensor, | ||
| + attention_mask: Tensor, | ||
| + key_value_states: Optional[Tensor] = None, | ||
| + inference_context: Optional[BaseInferenceContext] = None, | ||
| + rotary_pos_emb: Optional[Union[Tensor, Tuple[Tensor, Tensor]]] = None, | ||
| + rotary_pos_cos: Optional[Tensor] = None, | ||
| + rotary_pos_sin: Optional[Tensor] = None, | ||
| + attention_bias: Optional[Tensor] = None, | ||
| + packed_seq_params: Optional[PackedSeqParams] = None, | ||
| + sequence_len_offset: Optional[int] = None, | ||
| + *, | ||
| + inference_params: Optional[BaseInferenceContext] = None, | ||
| + ) -> Tuple[Tensor, Tensor]: | ||
| + """ | ||
| + Perform a forward pass through the attention module. | ||
| + | ||
| + Args: | ||
| + hidden_states (Tensor): Hidden states. | ||
| + attention_mask (Tensor): Attention mask. | ||
| + key_value_states (Optional[Tensor]): Key/value states (for cross attention). | ||
| + inference_context (Optional[BaseInferenceContext]): Inference context that manages | ||
| + KV cache. | ||
| + rotary_pos_emb (Optional[Union[Tensor, Tuple[Tensor, Tensor]]]): Rotary | ||
| + embedding tensor(s). | ||
| + rotary_pos_cos (Optional[Tensor]): Rotary embedding cosine. | ||
| + rotary_pos_sin (Optional[Tensor]): Rotary embedding sine. | ||
| + attention_bias (Optional[Tensor]): Attention bias. | ||
| + packed_seq_params (Optional[PackedSeqparams]): Parameters used for THD format. | ||
| + sequence_len_offset (Optional[int]): Sequence length offset used for | ||
| + inference CUDA graphs. | ||
| + | ||
| + Return: | ||
| + (Tuple[Tensor, Tensor]) Attention output and bias. | ||
| + | ||
| + """ | ||
| + # Check if we need to skip RoPE | ||
| + # no_rope is 0-indexed array and self.layer_number is 1-indexed | ||
| + assert ( | ||
| + attention_mask is not None | ||
| + ), "Flex attention is used for customed attention mask, which must be provided" | ||
| + | ||
| + no_rope = ( | ||
| + self.config.no_rope_freq[self.layer_number - 1] if self.config.no_rope_freq else False | ||
| + ) | ||
| + if no_rope: | ||
| + rotary_pos_emb = None | ||
| + | ||
| + inference_context = deprecate_inference_params(inference_context, inference_params) | ||
| + | ||
| + if inference_context and inference_context.is_dynamic_batching(): | ||
| + assert HAVE_FA3 or is_fa_min_version( | ||
| + "2.7.3" | ||
| + ), "flash attn verion v2.7.3 and above is required for dynamic batching." | ||
| + | ||
| + # hidden_states: [sq, b, h] | ||
| + if self.config.flash_decode and not self.training and inference_context is not None: | ||
| + rotary_pos_emb = None | ||
| + else: | ||
| + assert rotary_pos_cos is None and rotary_pos_sin is None | ||
| + | ||
| + # For self attention we just duplicate the rotary_pos_emb if it isn't already | ||
| + if rotary_pos_emb is not None and not isinstance(rotary_pos_emb, tuple): | ||
| + rotary_pos_emb = (rotary_pos_emb,) * 2 | ||
| + | ||
| + # ===================== | ||
| + # Query, Key, and Value | ||
| + # ===================== | ||
| + # Get the query, key and value tensors based on the type of attention - | ||
| + # self or cross attn. | ||
| + nvtx_range_push(suffix="qkv") | ||
| + query, key, value = self.get_query_key_value_tensors(hidden_states, key_value_states) | ||
| + nvtx_range_pop(suffix="qkv") | ||
| + | ||
| + | ||
| + # ================================================ | ||
| + # relative positional embedding (rotary embedding) | ||
| + # ================================================ | ||
| + nvtx_range_push(suffix="rotary_pos_emb") | ||
| + if rotary_pos_emb is not None and not self.config.flash_decode: | ||
| + q_pos_emb, k_pos_emb = rotary_pos_emb | ||
| + | ||
| + if packed_seq_params is not None: | ||
| + if packed_seq_params.cu_seqlens_q_padded is not None: | ||
| + cu_seqlens_q = packed_seq_params.cu_seqlens_q_padded | ||
| + else: | ||
| + cu_seqlens_q = packed_seq_params.cu_seqlens_q | ||
| + if packed_seq_params.cu_seqlens_kv_padded is not None: | ||
| + cu_seqlens_kv = packed_seq_params.cu_seqlens_kv_padded | ||
| + else: | ||
| + cu_seqlens_kv = packed_seq_params.cu_seqlens_kv | ||
| + else: | ||
| + cu_seqlens_q = cu_seqlens_kv = None | ||
| + | ||
| + if q_pos_emb is not None: | ||
| + # TODO VIJAY: simplify | ||
| + if inference_context is None or inference_context.is_static_batching(): | ||
| + query = apply_rotary_pos_emb( | ||
| + query, | ||
| + q_pos_emb, | ||
| + config=self.config, | ||
| + cu_seqlens=cu_seqlens_q, | ||
| + cp_group=self.pg_collection.cp, | ||
| + ) | ||
| + else: | ||
| + query = inference_context.apply_rotary_emb_query( | ||
| + query, q_pos_emb, self.config, cu_seqlens_q, self.pg_collection.cp | ||
| + ) | ||
| + if k_pos_emb is not None: | ||
| + key = apply_rotary_pos_emb( | ||
| + key, | ||
| + k_pos_emb, | ||
| + config=self.config, | ||
| + cu_seqlens=cu_seqlens_kv, | ||
| + cp_group=self.pg_collection.cp, | ||
| + ) | ||
| + | ||
| + # TODO, can apply positional embedding to value_layer so it has | ||
| + # absolute positional embedding. | ||
| + # otherwise, only relative positional embedding takes effect | ||
| + # value_layer = apply_rotary_pos_emb(value_layer, k_pos_emb) | ||
| + nvtx_range_pop(suffix="rotary_pos_emb") | ||
| + | ||
| + # ================================== | ||
| + # flex attention computation | ||
| + # ================================== | ||
| + query = query.permute( | ||
| + 1, 2, 0, 3 | ||
| + ).contiguous() # [seq_len, batch_size, num_heads/tp, head_dim] -> [batch_size, num_heads/tp, seq_len, head_dim] | ||
| + key = key.permute(1, 2, 0, 3).contiguous() | ||
| + value = value.permute(1, 2, 0, 3).contiguous() | ||
| + attention_mask = attention_mask.contiguous() | ||
| + core_attn_out, _ = ring_attn( | ||
| + query, key, value, attention_mask, group=get_context_parallel_group(), | ||
| + ) | ||
| + core_attn_out = core_attn_out.permute( | ||
| + 2, 0, 1, 3 | ||
| + ).contiguous() | ||
| + core_attn_out = core_attn_out.view(core_attn_out.shape[0], core_attn_out.shape[1], -1) | ||
| + | ||
| + | ||
| + # ================= | ||
| + # Output. [sq, b, h] | ||
| + # ================= | ||
| + | ||
| + output, bias = self.linear_proj(core_attn_out) | ||
| + | ||
| + return output, bias | ||
Oops, something went wrong.
Oops, something went wrong.
Add this suggestion to a batch that can be applied as a single commit.
This suggestion is invalid because no changes were made to the code.
Suggestions cannot be applied while the pull request is closed.
Suggestions cannot be applied while viewing a subset of changes.
Only one suggestion per line can be applied in a batch.
Add this suggestion to a batch that can be applied as a single commit.
Applying suggestions on deleted lines is not supported.
You must change the existing code in this line in order to create a valid suggestion.
Outdated suggestions cannot be applied.
This suggestion has been applied or marked resolved.
Suggestions cannot be applied from pending reviews.
Suggestions cannot be applied on multi-line comments.
Suggestions cannot be applied while the pull request is queued to merge.
Suggestion cannot be applied right now. Please check back later.
There was a problem hiding this comment.
Choose a reason for hiding this comment
The reason will be displayed to describe this comment to others. Learn more.
The variable
HAVE_FA3is used here but it's not defined or imported, which will lead to aNameErrorat runtime. It seems you intended to check for the availability of flash-attention v3.You should define
HAVE_FA3at the top of this file, similar to how it's done inmegatron/core/transformer/attention.py, by adding the followingtry-exceptblock: