Skip to content

[cuda] fix triu/tril int32 overflow for large matrices#164705

Closed
Aminsed wants to merge 2 commits intopytorch:mainfrom
Aminsed:fix/cuda-triu-tril-64bit-indexing
Closed

[cuda] fix triu/tril int32 overflow for large matrices#164705
Aminsed wants to merge 2 commits intopytorch:mainfrom
Aminsed:fix/cuda-triu-tril-64bit-indexing

Conversation

@Aminsed
Copy link
Contributor

@Aminsed Aminsed commented Oct 5, 2025

Fixes #136611

Cast blockIdx.x to int64_t before multiplication to prevent overflow when computing linear_idx for matrices larger than 2^31 elements.

cc @ngimel @ptrblck @eqy

Fixes pytorch#136611

Cast blockIdx.x to int64_t before multiplication to prevent overflow
when computing linear_idx for matrices larger than 2^31 elements.
@pytorch-bot
Copy link

pytorch-bot bot commented Oct 5, 2025

🔗 Helpful Links

🧪 See artifacts and rendered test results at hud.pytorch.org/pr/164705

Note: Links to docs will display an error until the docs builds have been completed.

❗ 1 Active SEVs

There are 1 currently active SEVs. If your PR is affected, please view them below:

✅ No Failures

As of commit d968f75 with merge base 321e602 (image):
💚 Looks good so far! There are no failures yet. 💚

This comment was automatically generated by Dr. CI and updates every 15 minutes.

@Aminsed
Copy link
Contributor Author

Aminsed commented Oct 5, 2025

@pytorchbot label "release notes: cuda"

@pytorch-bot pytorch-bot bot added the release notes: cuda release notes category label Oct 5, 2025
@eqy eqy added the ciflow/trunk Trigger trunk jobs on your pull request label Oct 5, 2025
@pytorch-bot
Copy link

pytorch-bot bot commented Oct 5, 2025

To add the ciflow label ciflow/trunk please first approve the workflows that are awaiting approval (scroll to the bottom of this page).

This helps ensure we don't trigger CI on this PR until it is actually authorized to do so. Please ping one of the reviewers if you do not have access to approve and run workflows.

@eqy eqy added the ciflow/h100 label Oct 5, 2025
@pytorch-bot pytorch-bot bot removed the ciflow/trunk Trigger trunk jobs on your pull request label Oct 5, 2025
@pytorch-bot
Copy link

pytorch-bot bot commented Oct 5, 2025

To add the ciflow label ciflow/h100 please first approve the workflows that are awaiting approval (scroll to the bottom of this page).

This helps ensure we don't trigger CI on this PR until it is actually authorized to do so. Please ping one of the reviewers if you do not have access to approve and run workflows.

@pytorch-bot pytorch-bot bot removed the ciflow/h100 label Oct 5, 2025
@eqy eqy added ciflow/trunk Trigger trunk jobs on your pull request ciflow/h100 labels Oct 5, 2025
@pytorch-bot
Copy link

pytorch-bot bot commented Oct 5, 2025

To add the ciflow label ciflow/trunk please first approve the workflows that are awaiting approval (scroll to the bottom of this page).

This helps ensure we don't trigger CI on this PR until it is actually authorized to do so. Please ping one of the reviewers if you do not have access to approve and run workflows.

@pytorch-bot
Copy link

pytorch-bot bot commented Oct 5, 2025

To add the ciflow label ciflow/h100 please first approve the workflows that are awaiting approval (scroll to the bottom of this page).

This helps ensure we don't trigger CI on this PR until it is actually authorized to do so. Please ping one of the reviewers if you do not have access to approve and run workflows.

@pytorch-bot pytorch-bot bot removed ciflow/trunk Trigger trunk jobs on your pull request ciflow/h100 labels Oct 5, 2025
@Aidyn-A Aidyn-A added ciflow/trunk Trigger trunk jobs on your pull request ciflow/h100 labels Oct 6, 2025
@ngimel
Copy link
Collaborator

ngimel commented Oct 6, 2025

Lint errors need to be fixed

@Aminsed
Copy link
Contributor Author

Aminsed commented Oct 19, 2025

@egy @ngimel Could you please merge as I'm not authorized.

@ngimel
Copy link
Collaborator

ngimel commented Oct 20, 2025

@pytorchbot merge

@pytorch-bot pytorch-bot bot added the ciflow/trunk Trigger trunk jobs on your pull request label Oct 20, 2025
@pytorchmergebot
Copy link
Collaborator

Merge started

Your change will be merged once all checks pass (ETA 0-4 Hours).

Learn more about merging in the wiki.

Questions? Feedback? Please reach out to the PyTorch DevX Team

Advanced Debugging
Check the merge workflow status
here

Comment on lines +9816 to +9824
@onlyCUDA
@largeTensorTest("40GB")
def test_triu_tril_large_matrix_64bit(self, device):
"""
Test triu/tril with large matrices requiring 64-bit indexing.
Regression test for https://github.com/pytorch/pytorch/issues/136611
"""
# 100k x 100k matrix with 10B elements requires 64-bit indexing
q_len = 100000
Copy link
Collaborator

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Copy link
Contributor Author

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

@onlyCUDA maybe?

Copy link
Collaborator

@nikitaved nikitaved Oct 20, 2025

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

Judging from the def, onlyCUDA does not seem to prevent ROCM, but there is skipCUDAIfRocm.

Copy link
Collaborator

Choose a reason for hiding this comment

The reason will be displayed to describe this comment to others. Learn more.

@skipCUDAIfROCm is what you're looking for. But I filed a DISABLED issue for this new test. We'll take a look as soon as possible as we continue to burn down skipped tests.

pytorchmergebot pushed a commit to renato-arantes/pytorch that referenced this pull request Oct 20, 2025
Fixes pytorch#136611

Cast blockIdx.x to int64_t before multiplication to prevent overflow when computing linear_idx for matrices larger than 2^31 elements.

Pull Request resolved: pytorch#164705
Approved by: https://github.com/eqy, https://github.com/ngimel
Chao1Han pushed a commit to Chao1Han/pytorch that referenced this pull request Oct 21, 2025
Fixes pytorch#136611

Cast blockIdx.x to int64_t before multiplication to prevent overflow when computing linear_idx for matrices larger than 2^31 elements.

Pull Request resolved: pytorch#164705
Approved by: https://github.com/eqy, https://github.com/ngimel
zhudada0120 pushed a commit to zhudada0120/pytorch that referenced this pull request Oct 22, 2025
Fixes pytorch#136611

Cast blockIdx.x to int64_t before multiplication to prevent overflow when computing linear_idx for matrices larger than 2^31 elements.

Pull Request resolved: pytorch#164705
Approved by: https://github.com/eqy, https://github.com/ngimel
Sign up for free to join this conversation on GitHub. Already have an account? Sign in to comment

Labels

ciflow/trunk Trigger trunk jobs on your pull request Merged open source release notes: cuda release notes category release notes: linalg_frontend release notes category

Projects

None yet

Development

Successfully merging this pull request may close these issues.

tensor.triu_(1) not working properly with large matrix

8 participants