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

optimize phi3 memory usage again #11848

Merged
merged 1 commit into from
Aug 19, 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
8 changes: 4 additions & 4 deletions python/llm/src/ipex_llm/transformers/models/phi3.py
Original file line number Diff line number Diff line change
Expand Up @@ -177,12 +177,12 @@ def attention_forward(
key_states = repeat_kv(key_states, self.num_key_value_groups)
value_states = repeat_kv(value_states, self.num_key_value_groups)

attn_weights = torch.matmul(query_states,
key_states.transpose(2, 3)) / math.sqrt(self.head_dim)
attn_weights = torch.matmul(query_states, key_states.transpose(2, 3))

# use inplaced div, add and softmax to avoid double attn_weights memory usage
attn_weights.div_(math.sqrt(self.head_dim))
if attention_mask is not None:
attn_weights = attn_weights + attention_mask

attn_weights.add_(attention_mask)
attn_weights = attention_softmax(attn_weights, self.training)

attn_weights = torch.nn.functional.dropout(attn_weights, p=self.attention_dropout,
Expand Down
Loading